Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Addition of Gradient Accumulation in LETR #15

Open
wants to merge 2 commits into
base: master
Choose a base branch
from
Open
Show file tree
Hide file tree
Changes from all commits
Commits
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
6 changes: 4 additions & 2 deletions src/args.py
Original file line number Diff line number Diff line change
Expand Up @@ -5,7 +5,9 @@ def get_args_parser():
parser = argparse.ArgumentParser('Set transformer detector', add_help=False)
parser.add_argument('--lr', default=1e-4, type=float)
parser.add_argument('--lr_backbone', default=1e-5, type=float)
parser.add_argument('--batch_size', default=2, type=int)
parser.add_argument('--batch_size', default=2, type=int)
parser.add_argument('--grad_accum_batches', default=1, type=int,
help="Number of batches to accumulate using gradient accumulation")
parser.add_argument('--weight_decay', default=1e-4, type=float)
parser.add_argument('--epochs', default=300, type=int)
parser.add_argument('--lr_drop', default=200, type=int)
Expand Down Expand Up @@ -115,4 +117,4 @@ def get_args_parser():
parser.add_argument('--eval', action='store_true')
parser.add_argument('--dataset', default='train', type=str, choices=('train', 'val'))

return parser
return parser
15 changes: 10 additions & 5 deletions src/engine.py
Original file line number Diff line number Diff line change
Expand Up @@ -25,6 +25,7 @@ def train_one_epoch(model, criterion, postprocessors, data_loader, optimizer, de


counter = 0
batch_idx = 0
torch.cuda.empty_cache()
for samples, targets in metric_logger.log_every(data_loader, print_freq, header):
samples = samples.to(device)
Expand Down Expand Up @@ -59,14 +60,18 @@ def train_one_epoch(model, criterion, postprocessors, data_loader, optimizer, de
print(loss_dict_reduced)
sys.exit(1)

optimizer.zero_grad()
losses /= args.grad_accum_batches
losses.backward()
if max_norm > 0:
torch.nn.utils.clip_grad_norm_(model.parameters(), max_norm)
optimizer.step()

if (batch_idx + 1) % args.grad_accum_batches == 0:
if args.clip_max_norm > 0:
torch.nn.utils.clip_grad_norm_(model.parameters(), args.clip_max_norm)
optimizer.step()
optimizer.zero_grad()

metric_logger.update(loss=loss_value, **loss_dict_reduced_scaled, **loss_dict_reduced_unscaled)
metric_logger.update(lr=optimizer.param_groups[0]["lr"])
batch_idx +=1
# gather the stats from all processes
metric_logger.synchronize_between_processes()
print("Averaged stats:", metric_logger)
Expand Down Expand Up @@ -163,4 +168,4 @@ def evaluate(model, criterion, postprocessors, data_loader, base_ds, device, out
# accumulate predictions from all images
stats = {k: meter.global_avg for k, meter in metric_logger.meters.items()}

return stats
return stats