You can not select more than 25 topics Topics must start with a letter or number, can include dashes ('-') and can be up to 35 characters long.

276 lines
10 KiB

  1. import os
  2. import time
  3. import argparse
  4. import math
  5. import torch
  6. from distributed import DistributedDataParallel
  7. from torch.utils.data.distributed import DistributedSampler
  8. from torch.nn import DataParallel
  9. from torch.utils.data import DataLoader
  10. from fp16_optimizer import FP16_Optimizer
  11. from model import Tacotron2
  12. from data_utils import TextMelLoader, TextMelCollate
  13. from loss_function import Tacotron2Loss
  14. from logger import Tacotron2Logger
  15. from hparams import create_hparams
  16. def batchnorm_to_float(module):
  17. """Converts batch norm modules to FP32"""
  18. if isinstance(module, torch.nn.modules.batchnorm._BatchNorm):
  19. module.float()
  20. for child in module.children():
  21. batchnorm_to_float(child)
  22. return module
  23. def reduce_tensor(tensor, num_gpus):
  24. rt = tensor.clone()
  25. torch.distributed.all_reduce(rt, op=torch.distributed.reduce_op.SUM)
  26. rt /= num_gpus
  27. return rt
  28. def init_distributed(hparams, n_gpus, rank, group_name):
  29. assert torch.cuda.is_available(), "Distributed mode requires CUDA."
  30. print("Initializing distributed")
  31. # Set cuda device so everything is done on the right GPU.
  32. torch.cuda.set_device(rank % torch.cuda.device_count())
  33. # Initialize distributed communication
  34. torch.distributed.init_process_group(
  35. backend=hparams.dist_backend, init_method=hparams.dist_url,
  36. world_size=n_gpus, rank=rank, group_name=group_name)
  37. print("Done initializing distributed")
  38. def prepare_dataloaders(hparams):
  39. # Get data, data loaders and collate function ready
  40. trainset = TextMelLoader(hparams.training_files, hparams)
  41. valset = TextMelLoader(hparams.validation_files, hparams)
  42. collate_fn = TextMelCollate(hparams.n_frames_per_step)
  43. train_sampler = DistributedSampler(trainset) \
  44. if hparams.distributed_run else None
  45. train_loader = DataLoader(trainset, num_workers=1, shuffle=False,
  46. sampler=train_sampler,
  47. batch_size=hparams.batch_size, pin_memory=False,
  48. drop_last=True, collate_fn=collate_fn)
  49. return train_loader, valset, collate_fn
  50. def prepare_directories_and_logger(output_directory, log_directory, rank):
  51. if rank == 0:
  52. if not os.path.isdir(output_directory):
  53. os.makedirs(output_directory)
  54. os.chmod(output_directory, 0o775)
  55. logger = Tacotron2Logger(os.path.join(output_directory, log_directory))
  56. else:
  57. logger = None
  58. return logger
  59. def load_model(hparams):
  60. model = Tacotron2(hparams).cuda()
  61. model = batchnorm_to_float(model.half()) if hparams.fp16_run else model
  62. if hparams.distributed_run:
  63. model = DistributedDataParallel(model)
  64. elif torch.cuda.device_count() > 1:
  65. model = DataParallel(model)
  66. return model
  67. def warm_start_model(checkpoint_path, model):
  68. assert os.path.isfile(checkpoint_path)
  69. print("Warm starting model from checkpoint '{}'".format(checkpoint_path))
  70. checkpoint_dict = torch.load(checkpoint_path, map_location='cpu')
  71. model.load_state_dict(checkpoint_dict['state_dict'])
  72. return model
  73. def load_checkpoint(checkpoint_path, model, optimizer):
  74. assert os.path.isfile(checkpoint_path)
  75. print("Loading checkpoint '{}'".format(checkpoint_path))
  76. checkpoint_dict = torch.load(checkpoint_path, map_location='cpu')
  77. model.load_state_dict(checkpoint_dict['state_dict'])
  78. optimizer.load_state_dict(checkpoint_dict['optimizer'])
  79. learning_rate = checkpoint_dict['learning_rate']
  80. iteration = checkpoint_dict['iteration']
  81. print("Loaded checkpoint '{}' from iteration {}" .format(
  82. checkpoint_path, iteration))
  83. return model, optimizer, learning_rate, iteration
  84. def save_checkpoint(model, optimizer, learning_rate, iteration, filepath):
  85. print("Saving model and optimizer state at iteration {} to {}".format(
  86. iteration, filepath))
  87. torch.save({'iteration': iteration,
  88. 'state_dict': model.state_dict(),
  89. 'optimizer': optimizer.state_dict(),
  90. 'learning_rate': learning_rate}, filepath)
  91. def validate(model, criterion, valset, iteration, batch_size, n_gpus,
  92. collate_fn, logger, distributed_run, rank):
  93. """Handles all the validation scoring and printing"""
  94. model.eval()
  95. with torch.no_grad():
  96. val_sampler = DistributedSampler(valset) if distributed_run else None
  97. val_loader = DataLoader(valset, sampler=val_sampler, num_workers=1,
  98. shuffle=False, batch_size=batch_size,
  99. pin_memory=False, collate_fn=collate_fn)
  100. val_loss = 0.0
  101. for i, batch in enumerate(val_loader):
  102. x, y = model.parse_batch(batch)
  103. y_pred = model(x)
  104. loss = criterion(y_pred, y)
  105. reduced_val_loss = reduce_tensor(loss.data, n_gpus)[0] \
  106. if distributed_run else loss.data[0]
  107. val_loss += reduced_val_loss
  108. val_loss = val_loss / (i + 1)
  109. model.train()
  110. return val_loss
  111. def train(output_directory, log_directory, checkpoint_path, warm_start, n_gpus,
  112. rank, group_name, hparams):
  113. """Training and validation logging results to tensorboard and stdout
  114. Params
  115. ------
  116. output_directory (string): directory to save checkpoints
  117. log_directory (string) directory to save tensorboard logs
  118. checkpoint_path(string): checkpoint path
  119. n_gpus (int): number of gpus
  120. rank (int): rank of current gpu
  121. hparams (object): comma separated list of "name=value" pairs.
  122. """
  123. if hparams.distributed_run:
  124. init_distributed(hparams, n_gpus, rank, group_name)
  125. torch.manual_seed(hparams.seed)
  126. torch.cuda.manual_seed(hparams.seed)
  127. model = load_model(hparams)
  128. learning_rate = hparams.learning_rate
  129. optimizer = torch.optim.Adam(model.parameters(), lr=learning_rate,
  130. weight_decay=hparams.weight_decay)
  131. if hparams.fp16_run:
  132. optimizer = FP16_Optimizer(
  133. optimizer, dynamic_loss_scale=hparams.dynamic_loss_scaling)
  134. criterion = Tacotron2Loss()
  135. logger = prepare_directories_and_logger(
  136. output_directory, log_directory, rank)
  137. train_loader, valset, collate_fn = prepare_dataloaders(hparams)
  138. # Load checkpoint if one exists
  139. iteration = 0
  140. epoch_offset = 0
  141. if checkpoint_path is not None:
  142. if warm_start:
  143. model = warm_start_model(checkpoint_path, model)
  144. else:
  145. model, optimizer, learning_rate, iteration = load_checkpoint(
  146. checkpoint_path, model, optimizer)
  147. iteration += 1 # next iteration is iteration + 1
  148. epoch_offset = max(0, int(iteration / len(train_loader)))
  149. model.train()
  150. # ================ MAIN TRAINNIG LOOP! ===================
  151. for epoch in range(epoch_offset, hparams.epochs):
  152. print("Epoch: {}".format(epoch))
  153. for i, batch in enumerate(train_loader):
  154. start = time.perf_counter()
  155. for param_group in optimizer.param_groups:
  156. param_group['lr'] = learning_rate
  157. model.zero_grad()
  158. x, y = model.parse_batch(batch)
  159. y_pred = model(x)
  160. loss = criterion(y_pred, y)
  161. reduced_loss = reduce_tensor(loss.data, n_gpus)[0] \
  162. if hparams.distributed_run else loss.data[0]
  163. if hparams.fp16_run:
  164. optimizer.backward(loss)
  165. grad_norm = optimizer.clip_fp32_grads(hparams.grad_clip_thresh)
  166. else:
  167. loss.backward()
  168. grad_norm = torch.nn.utils.clip_grad_norm(
  169. model.parameters(), hparams.grad_clip_thresh)
  170. optimizer.step()
  171. overflow = optimizer.overflow if hparams.fp16_run else False
  172. if not overflow and not math.isnan(reduced_loss) and rank == 0:
  173. duration = time.perf_counter() - start
  174. print("Train loss {} {:.6f} Grad Norm {:.6f} {:.2f}s/it".format(
  175. iteration, reduced_loss, grad_norm, duration))
  176. logger.log_training(
  177. reduced_loss, grad_norm, learning_rate, duration, iteration)
  178. if not overflow and (iteration % hparams.iters_per_checkpoint == 0):
  179. reduced_val_loss = validate(
  180. model, criterion, valset, iteration, hparams.batch_size,
  181. n_gpus, collate_fn, logger, hparams.distributed_run, rank)
  182. if rank == 0:
  183. print("Validation loss {}: {:9f} ".format(
  184. iteration, reduced_val_loss))
  185. logger.log_validation(
  186. reduced_val_loss, model, y, y_pred, iteration)
  187. checkpoint_path = os.path.join(
  188. output_directory, "checkpoint_{}".format(iteration))
  189. save_checkpoint(model, optimizer, learning_rate, iteration,
  190. checkpoint_path)
  191. iteration += 1
  192. if __name__ == '__main__':
  193. parser = argparse.ArgumentParser()
  194. parser.add_argument('-o', '--output_directory', type=str,
  195. help='directory to save checkpoints')
  196. parser.add_argument('-l', '--log_directory', type=str,
  197. help='directory to save tensorboard logs')
  198. parser.add_argument('-c', '--checkpoint_path', type=str, default=None,
  199. required=False, help='checkpoint path')
  200. parser.add_argument('--warm_start', action='store_true',
  201. help='load the model only (warm start)')
  202. parser.add_argument('--n_gpus', type=int, default=1,
  203. required=False, help='number of gpus')
  204. parser.add_argument('--rank', type=int, default=0,
  205. required=False, help='rank of current gpu')
  206. parser.add_argument('--group_name', type=str, default='group_name',
  207. required=False, help='Distributed group name')
  208. parser.add_argument('--hparams', type=str,
  209. required=False, help='comma separated name=value pairs')
  210. args = parser.parse_args()
  211. hparams = create_hparams(args.hparams)
  212. torch.backends.cudnn.enabled = hparams.cudnn_enabled
  213. torch.backends.cudnn.benchmark = hparams.cudnn_benchmark
  214. print("FP16 Run:", hparams.fp16_run)
  215. print("Dynamic Loss Scaling", hparams.dynamic_loss_scaling)
  216. print("Distributed Run:", hparams.distributed_run)
  217. print("cuDNN Enabled:", hparams.cudnn_enabled)
  218. print("cuDNN Benchmark:", hparams.cudnn_benchmark)
  219. train(args.output_directory, args.log_directory, args.checkpoint_path,
  220. args.warm_start, args.n_gpus, args.rank, args.group_name, hparams)