engine.py 22 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557
  1. import torch
  2. import torch.distributed as dist
  3. import os
  4. import random
  5. # ----------------- Extra Components -----------------
  6. from utils import distributed_utils
  7. from utils.misc import MetricLogger, SmoothedValue
  8. from utils.vis_tools import vis_data
  9. # ----------------- Optimizer & LrScheduler Components -----------------
  10. from utils.solver.optimizer import build_yolo_optimizer, build_rtdetr_optimizer
  11. from utils.solver.lr_scheduler import LinearWarmUpLrScheduler, build_lr_scheduler
  12. class YoloTrainer(object):
  13. def __init__(self,
  14. # Basic parameters
  15. args,
  16. cfg,
  17. device,
  18. # Model parameters
  19. model,
  20. model_ema,
  21. criterion,
  22. # Data parameters
  23. train_transform,
  24. val_transform,
  25. dataset,
  26. train_loader,
  27. evaluator,
  28. ):
  29. # ------------------- basic parameters -------------------
  30. self.args = args
  31. self.cfg = cfg
  32. self.epoch = 0
  33. self.best_map = -1.
  34. self.device = device
  35. self.criterion = criterion
  36. self.heavy_eval = False
  37. self.model_ema = model_ema
  38. # weak augmentatino stage
  39. self.second_stage = False
  40. self.second_stage_epoch = cfg.no_aug_epoch
  41. # path to save model
  42. self.path_to_save = os.path.join(args.save_folder, args.dataset, args.model)
  43. os.makedirs(self.path_to_save, exist_ok=True)
  44. # ---------------------------- Transform ----------------------------
  45. self.train_transform = train_transform
  46. self.val_transform = val_transform
  47. # ---------------------------- Dataset & Dataloader ----------------------------
  48. self.dataset = dataset
  49. self.train_loader = train_loader
  50. # ---------------------------- Evaluator ----------------------------
  51. self.evaluator = evaluator
  52. # ---------------------------- Build Grad. Scaler ----------------------------
  53. self.scaler = torch.cuda.amp.GradScaler(enabled=args.fp16)
  54. # ---------------------------- Build Optimizer ----------------------------
  55. self.grad_accumulate = max(64 // args.batch_size, 1)
  56. cfg.base_lr = cfg.per_image_lr * args.batch_size * self.grad_accumulate
  57. cfg.min_lr = cfg.base_lr * cfg.min_lr_ratio
  58. self.optimizer, self.start_epoch = build_yolo_optimizer(cfg, model, args.resume)
  59. # ---------------------------- Build LR Scheduler ----------------------------
  60. warmup_iters = cfg.warmup_epoch * len(self.train_loader)
  61. self.lr_scheduler_warmup = LinearWarmUpLrScheduler(warmup_iters, cfg.base_lr, cfg.warmup_bias_lr)
  62. self.lr_scheduler = build_lr_scheduler(cfg, self.optimizer, args.resume)
  63. self.best_map = cfg.best_map
  64. def train(self, model):
  65. for epoch in range(self.start_epoch, self.cfg.max_epoch):
  66. if self.args.distributed:
  67. self.train_loader.batch_sampler.sampler.set_epoch(epoch)
  68. # check second stage
  69. if epoch >= (self.cfg.max_epoch - self.second_stage_epoch - 1) and not self.second_stage:
  70. self.check_second_stage()
  71. # save model of the last mosaic epoch
  72. weight_name = '{}_last_mosaic_epoch.pth'.format(self.args.model)
  73. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  74. print('Saving state of the last Mosaic epoch-{}.'.format(self.epoch))
  75. torch.save({'model': model.state_dict(),
  76. 'mAP': round(self.evaluator.map*100, 1),
  77. 'optimizer': self.optimizer.state_dict(),
  78. 'epoch': self.epoch,
  79. 'args': self.args},
  80. checkpoint_path)
  81. # train one epoch
  82. self.epoch = epoch
  83. self.train_one_epoch(model)
  84. # LR Schedule
  85. if (epoch + 1) > self.cfg.warmup_epoch:
  86. self.lr_scheduler.step()
  87. # eval one epoch
  88. if self.heavy_eval:
  89. model_eval = model.module if self.args.distributed else model
  90. self.eval(model_eval)
  91. else:
  92. model_eval = model.module if self.args.distributed else model
  93. if (epoch % self.cfg.eval_epoch) == 0 or (epoch == self.cfg.max_epoch - 1):
  94. self.eval(model_eval)
  95. if self.args.debug:
  96. print("For debug mode, we only train 1 epoch")
  97. break
  98. def eval(self, model):
  99. # set eval mode
  100. model.eval()
  101. model_eval = model if self.model_ema is None else self.model_ema.ema
  102. cur_map = -1.
  103. to_save = False
  104. if distributed_utils.is_main_process():
  105. if self.evaluator is None:
  106. print('No evaluator ... save model and go on training.')
  107. to_save = True
  108. weight_name = '{}_no_eval.pth'.format(self.args.model)
  109. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  110. else:
  111. print('Eval ...')
  112. # Evaluate
  113. with torch.no_grad():
  114. self.evaluator.evaluate(model_eval)
  115. cur_map = self.evaluator.map
  116. if cur_map > self.best_map:
  117. # update best-map
  118. self.best_map = cur_map
  119. to_save = True
  120. # Save model
  121. if to_save:
  122. print('Saving state, epoch:', self.epoch)
  123. weight_name = '{}_best.pth'.format(self.args.model)
  124. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  125. state_dicts = {
  126. 'model': model_eval.state_dict(),
  127. 'mAP': round(cur_map*100, 1),
  128. 'optimizer': self.optimizer.state_dict(),
  129. 'lr_scheduler': self.lr_scheduler.state_dict(),
  130. 'epoch': self.epoch,
  131. 'args': self.args,
  132. }
  133. if self.model_ema is not None:
  134. state_dicts["ema_updates"] = self.model_ema.updates
  135. torch.save(state_dicts, checkpoint_path)
  136. if self.args.distributed:
  137. # wait for all processes to synchronize
  138. dist.barrier()
  139. # set train mode.
  140. model.train()
  141. def train_one_epoch(self, model):
  142. metric_logger = MetricLogger(delimiter=" ")
  143. metric_logger.add_meter('lr', SmoothedValue(window_size=1, fmt='{value:.6f}'))
  144. metric_logger.add_meter('size', SmoothedValue(window_size=1, fmt='{value:d}'))
  145. metric_logger.add_meter('gnorm', SmoothedValue(window_size=1, fmt='{value:.1f}'))
  146. header = 'Epoch: [{} / {}]'.format(self.epoch, self.cfg.max_epoch)
  147. epoch_size = len(self.train_loader)
  148. print_freq = 10
  149. gnorm = 0.0
  150. # basic parameters
  151. epoch_size = len(self.train_loader)
  152. img_size = self.cfg.train_img_size
  153. nw = epoch_size * self.cfg.warmup_epoch
  154. # Train one epoch
  155. for iter_i, (images, targets) in enumerate(metric_logger.log_every(self.train_loader, print_freq, header)):
  156. ni = iter_i + self.epoch * epoch_size
  157. # Warmup
  158. if nw > 0 and ni < nw:
  159. self.lr_scheduler_warmup(ni, self.optimizer)
  160. elif ni == nw:
  161. print("Warmup stage is over.")
  162. self.lr_scheduler_warmup.set_lr(self.optimizer, self.cfg.base_lr)
  163. # To device
  164. images = images.to(self.device, non_blocking=True).float()
  165. # Multi scale
  166. images, targets, img_size = self.rescale_image_targets(
  167. images, targets, self.cfg.max_stride, self.cfg.multi_scale)
  168. # Visualize train targets
  169. if self.args.vis_tgt:
  170. vis_data(images,
  171. targets,
  172. self.cfg.num_classes,
  173. self.cfg.normalize_coords,
  174. self.train_transform.color_format,
  175. self.cfg.pixel_mean,
  176. self.cfg.pixel_std,
  177. self.cfg.box_format)
  178. # Inference
  179. with torch.cuda.amp.autocast(enabled=self.args.fp16):
  180. outputs = model(images)
  181. # Compute loss
  182. loss_dict = self.criterion(outputs=outputs, targets=targets)
  183. losses = loss_dict['losses']
  184. losses /= self.grad_accumulate
  185. loss_dict_reduced = distributed_utils.reduce_dict(loss_dict)
  186. # Backward
  187. self.scaler.scale(losses).backward()
  188. # Optimize
  189. if (iter_i + 1) % self.grad_accumulate == 0:
  190. if self.cfg.clip_max_norm > 0:
  191. self.scaler.unscale_(self.optimizer)
  192. gnorm = torch.nn.utils.clip_grad_norm_(model.parameters(), max_norm=self.cfg.clip_max_norm)
  193. self.scaler.step(self.optimizer)
  194. self.scaler.update()
  195. self.optimizer.zero_grad()
  196. # ModelEMA
  197. if self.model_ema is not None:
  198. self.model_ema.update(model)
  199. # Update log
  200. metric_logger.update(**loss_dict_reduced)
  201. metric_logger.update(lr=self.optimizer.param_groups[2]["lr"])
  202. metric_logger.update(size=img_size)
  203. metric_logger.update(gnorm=gnorm)
  204. if self.args.debug:
  205. print("For debug mode, we only train 1 iteration")
  206. break
  207. # Gather the stats from all processes
  208. metric_logger.synchronize_between_processes()
  209. print("Averaged stats:", metric_logger)
  210. def rescale_image_targets(self, images, targets, max_stride, multi_scale_range=[0.5, 1.5]):
  211. """
  212. Deployed for Multi scale trick.
  213. """
  214. # During training phase, the shape of input image is square.
  215. old_img_size = images.shape[-1]
  216. min_img_size = old_img_size * multi_scale_range[0]
  217. max_img_size = old_img_size * multi_scale_range[1]
  218. # Choose a new image size
  219. new_img_size = random.randrange(min_img_size, max_img_size + max_stride, max_stride)
  220. # Resize
  221. if new_img_size != old_img_size:
  222. # interpolate
  223. images = torch.nn.functional.interpolate(
  224. input=images,
  225. size=new_img_size,
  226. mode='bilinear',
  227. align_corners=False)
  228. # rescale targets
  229. if not self.cfg.normalize_coords:
  230. for tgt in targets:
  231. boxes = tgt["boxes"].clone()
  232. labels = tgt["labels"].clone()
  233. boxes = torch.clamp(boxes, 0, old_img_size)
  234. # rescale box
  235. boxes[:, [0, 2]] = boxes[:, [0, 2]] / old_img_size * new_img_size
  236. boxes[:, [1, 3]] = boxes[:, [1, 3]] / old_img_size * new_img_size
  237. # refine tgt
  238. tgt_boxes_wh = boxes[..., 2:] - boxes[..., :2]
  239. min_tgt_size = torch.min(tgt_boxes_wh, dim=-1)[0]
  240. keep = (min_tgt_size >= 8)
  241. tgt["boxes"] = boxes[keep]
  242. tgt["labels"] = labels[keep]
  243. return images, targets, new_img_size
  244. def check_second_stage(self):
  245. # set second stage
  246. print('============== Second stage of Training ==============')
  247. self.second_stage = True
  248. self.heavy_eval = True
  249. # close mosaic augmentation
  250. if self.train_loader.dataset.mosaic_prob > 0.:
  251. print(' - Close < Mosaic Augmentation > ...')
  252. self.train_loader.dataset.mosaic_prob = 0.
  253. # close mixup augmentation
  254. if self.train_loader.dataset.mixup_prob > 0.:
  255. print(' - Close < Mixup Augmentation > ...')
  256. self.train_loader.dataset.mixup_prob = 0.
  257. # close copy-paste augmentation
  258. if self.train_loader.dataset.copy_paste > 0.:
  259. print(' - Close < Copy-paste Augmentation > ...')
  260. self.train_loader.dataset.copy_paste = 0.
  261. class RTDetrTrainer(object):
  262. def __init__(self,
  263. # Basic parameters
  264. args,
  265. cfg,
  266. device,
  267. # Model parameters
  268. model,
  269. model_ema,
  270. criterion,
  271. # Data parameters
  272. train_transform,
  273. val_transform,
  274. dataset,
  275. train_loader,
  276. evaluator,
  277. ):
  278. # ------------------- basic parameters -------------------
  279. self.args = args
  280. self.cfg = cfg
  281. self.epoch = 0
  282. self.best_map = -1.
  283. self.device = device
  284. self.criterion = criterion
  285. self.heavy_eval = False
  286. self.model_ema = model_ema
  287. # path to save model
  288. self.path_to_save = os.path.join(args.save_folder, args.dataset, args.model)
  289. os.makedirs(self.path_to_save, exist_ok=True)
  290. # ---------------------------- Transform ----------------------------
  291. self.train_transform = train_transform
  292. self.val_transform = val_transform
  293. # ---------------------------- Dataset & Dataloader ----------------------------
  294. self.dataset = dataset
  295. self.train_loader = train_loader
  296. # ---------------------------- Evaluator ----------------------------
  297. self.evaluator = evaluator
  298. # ---------------------------- Build Grad. Scaler ----------------------------
  299. self.scaler = torch.cuda.amp.GradScaler(enabled=args.fp16)
  300. # ---------------------------- Build Optimizer ----------------------------
  301. self.grad_accumulate = max(16 // args.batch_size, 1)
  302. cfg.base_lr = cfg.per_image_lr * args.batch_size * self.grad_accumulate
  303. cfg.min_lr = cfg.base_lr * cfg.min_lr_ratio
  304. self.optimizer, self.start_epoch = build_rtdetr_optimizer(cfg, model, args.resume)
  305. # ---------------------------- Build LR Scheduler ----------------------------
  306. self.wp_lr_scheduler = LinearWarmUpLrScheduler(cfg.warmup_iters, cfg.base_lr)
  307. self.lr_scheduler = build_lr_scheduler(cfg, self.optimizer, args.resume)
  308. def train(self, model):
  309. for epoch in range(self.start_epoch, self.cfg.max_epoch):
  310. if self.args.distributed:
  311. self.train_loader.batch_sampler.sampler.set_epoch(epoch)
  312. # train one epoch
  313. self.epoch = epoch
  314. self.train_one_epoch(model)
  315. # LR Scheduler
  316. self.lr_scheduler.step()
  317. # eval one epoch
  318. if self.heavy_eval:
  319. model_eval = model.module if self.args.distributed else model
  320. self.eval(model_eval)
  321. else:
  322. model_eval = model.module if self.args.distributed else model
  323. if (epoch % self.cfg.eval_epoch) == 0 or (epoch == self.cfg.max_epoch - 1):
  324. self.eval(model_eval)
  325. if self.args.debug:
  326. print("For debug mode, we only train 1 epoch")
  327. break
  328. def eval(self, model):
  329. # set eval mode
  330. model.eval()
  331. model_eval = model if self.model_ema is None else self.model_ema.ema
  332. cur_map = -1.
  333. to_save = False
  334. if distributed_utils.is_main_process():
  335. if self.evaluator is None:
  336. print('No evaluator ... save model and go on training.')
  337. to_save = True
  338. weight_name = '{}_no_eval.pth'.format(self.args.model)
  339. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  340. else:
  341. print('Eval ...')
  342. # Evaluate
  343. with torch.no_grad():
  344. self.evaluator.evaluate(model_eval)
  345. cur_map = self.evaluator.map
  346. if cur_map > self.best_map:
  347. # update best-map
  348. self.best_map = cur_map
  349. to_save = True
  350. # Save model
  351. if to_save:
  352. print('Saving state, epoch:', self.epoch)
  353. weight_name = '{}_best.pth'.format(self.args.model)
  354. checkpoint_path = os.path.join(self.path_to_save, weight_name)
  355. state_dicts = {
  356. 'model': model_eval.state_dict(),
  357. 'mAP': round(cur_map*100, 1),
  358. 'optimizer': self.optimizer.state_dict(),
  359. 'lr_scheduler': self.lr_scheduler.state_dict(),
  360. 'epoch': self.epoch,
  361. 'args': self.args,
  362. }
  363. if self.model_ema is not None:
  364. state_dicts["ema_updates"] = self.model_ema.updates
  365. torch.save(state_dicts, checkpoint_path)
  366. if self.args.distributed:
  367. # wait for all processes to synchronize
  368. dist.barrier()
  369. # set train mode.
  370. model.train()
  371. def train_one_epoch(self, model):
  372. metric_logger = MetricLogger(delimiter=" ")
  373. metric_logger.add_meter('lr', SmoothedValue(window_size=1, fmt='{value:.6f}'))
  374. metric_logger.add_meter('size', SmoothedValue(window_size=1, fmt='{value:d}'))
  375. metric_logger.add_meter('gnorm', SmoothedValue(window_size=1, fmt='{value:.1f}'))
  376. header = 'Epoch: [{} / {}]'.format(self.epoch, self.cfg.max_epoch)
  377. epoch_size = len(self.train_loader)
  378. print_freq = 10
  379. gnorm = 0.0
  380. # basic parameters
  381. epoch_size = len(self.train_loader)
  382. img_size = self.cfg.train_img_size
  383. nw = self.cfg.warmup_iters
  384. lr_warmup_stage = True
  385. # Train one epoch
  386. for iter_i, (images, targets) in enumerate(metric_logger.log_every(self.train_loader, print_freq, header)):
  387. ni = iter_i + self.epoch * epoch_size
  388. # WarmUp
  389. if lr_warmup_stage:
  390. if ni % self.grad_accumulate == 0:
  391. ni = ni // self.grad_accumulate
  392. if ni < nw:
  393. self.wp_lr_scheduler(ni, self.optimizer)
  394. elif ni == nw and lr_warmup_stage:
  395. print('Warmup stage is over.')
  396. lr_warmup_stage = False
  397. self.wp_lr_scheduler.set_lr(self.optimizer, self.cfg.base_lr)
  398. # To device
  399. images = images.to(self.device, non_blocking=True).float()
  400. for tgt in targets:
  401. tgt['boxes'] = tgt['boxes'].to(self.device)
  402. tgt['labels'] = tgt['labels'].to(self.device)
  403. # Multi scale
  404. images, targets, img_size = self.rescale_image_targets(
  405. images, targets, self.cfg.max_stride, self.cfg.multi_scale)
  406. # Visualize train targets
  407. if self.args.vis_tgt:
  408. vis_data(images,
  409. targets,
  410. self.cfg.num_classes,
  411. self.cfg.normalize_coords,
  412. self.train_transform.color_format,
  413. self.cfg.pixel_mean,
  414. self.cfg.pixel_std,
  415. self.cfg.box_format)
  416. # Inference
  417. with torch.cuda.amp.autocast(enabled=self.args.fp16):
  418. outputs = model(images, targets)
  419. # Compute loss
  420. loss_dict = self.criterion(outputs, targets)
  421. losses = sum(loss_dict.values())
  422. losses /= self.grad_accumulate
  423. loss_dict_reduced = distributed_utils.reduce_dict(loss_dict)
  424. # Backward
  425. self.scaler.scale(losses).backward()
  426. # Optimize
  427. if (iter_i + 1) % self.grad_accumulate == 0:
  428. if self.cfg.clip_max_norm > 0:
  429. self.scaler.unscale_(self.optimizer)
  430. gnorm = torch.nn.utils.clip_grad_norm_(model.parameters(), max_norm=self.cfg.clip_max_norm)
  431. self.scaler.step(self.optimizer)
  432. self.scaler.update()
  433. self.optimizer.zero_grad()
  434. # ModelEMA
  435. if self.model_ema is not None:
  436. self.model_ema.update(model)
  437. # Update log
  438. metric_logger.update(**loss_dict_reduced)
  439. metric_logger.update(lr=self.optimizer.param_groups[2]["lr"])
  440. metric_logger.update(size=img_size)
  441. metric_logger.update(gnorm=gnorm)
  442. if self.args.debug:
  443. print("For debug mode, we only train 1 iteration")
  444. break
  445. def rescale_image_targets(self, images, targets, max_stride, multi_scale_range=[0.5, 1.5]):
  446. """
  447. Deployed for Multi scale trick.
  448. """
  449. # During training phase, the shape of input image is square.
  450. old_img_size = images.shape[-1]
  451. min_img_size = old_img_size * multi_scale_range[0]
  452. max_img_size = old_img_size * multi_scale_range[1]
  453. # Choose a new image size
  454. new_img_size = random.randrange(min_img_size, max_img_size + max_stride, max_stride)
  455. # Resize
  456. if new_img_size != old_img_size:
  457. # interpolate
  458. images = torch.nn.functional.interpolate(
  459. input=images,
  460. size=new_img_size,
  461. mode='bilinear',
  462. align_corners=False)
  463. return images, targets, new_img_size
  464. # Build Trainer
  465. def build_trainer(args, cfg, device, model, model_ema, criterion, train_transform, val_transform, dataset, train_loader, evaluator):
  466. # ----------------------- Det trainers -----------------------
  467. if cfg.trainer == 'yolo':
  468. return YoloTrainer(args, cfg, device, model, model_ema, criterion, train_transform, val_transform, dataset, train_loader, evaluator)
  469. elif cfg.trainer == 'rtdetr':
  470. return RTDetrTrainer(args, cfg, device, model, model_ema, criterion, train_transform, val_transform, dataset, train_loader, evaluator)
  471. else:
  472. raise NotImplementedError(cfg.trainer)