1
0

flow.py 6.3 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149
  1. # Copyright (c) 2024 Alibaba Inc (authors: Xiang Lyu, Zhihao Du)
  2. #
  3. # Licensed under the Apache License, Version 2.0 (the "License");
  4. # you may not use this file except in compliance with the License.
  5. # You may obtain a copy of the License at
  6. #
  7. # http://www.apache.org/licenses/LICENSE-2.0
  8. #
  9. # Unless required by applicable law or agreed to in writing, software
  10. # distributed under the License is distributed on an "AS IS" BASIS,
  11. # WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
  12. # See the License for the specific language governing permissions and
  13. # limitations under the License.
  14. import logging
  15. import random
  16. from typing import Dict, Optional
  17. import torch
  18. import torch.nn as nn
  19. from torch.nn import functional as F
  20. from omegaconf import DictConfig
  21. from cosyvoice.utils.mask import make_pad_mask
  22. class MaskedDiffWithXvec(torch.nn.Module):
  23. def __init__(self,
  24. input_size: int = 512,
  25. output_size: int = 80,
  26. spk_embed_dim: int = 192,
  27. output_type: str = "mel",
  28. vocab_size: int = 4096,
  29. input_frame_rate: int = 50,
  30. only_mask_loss: bool = True,
  31. encoder: torch.nn.Module = None,
  32. length_regulator: torch.nn.Module = None,
  33. decoder: torch.nn.Module = None,
  34. decoder_conf: Dict = {'in_channels': 240, 'out_channel': 80, 'spk_emb_dim': 80, 'n_spks': 1,
  35. 'cfm_params': DictConfig({'sigma_min': 1e-06, 'solver': 'euler', 't_scheduler': 'cosine',
  36. 'training_cfg_rate': 0.2, 'inference_cfg_rate': 0.7, 'reg_loss_type': 'l1'}),
  37. 'decoder_params': {'channels': [256, 256], 'dropout': 0.0, 'attention_head_dim': 64,
  38. 'n_blocks': 4, 'num_mid_blocks': 12, 'num_heads': 8, 'act_fn': 'gelu'}},
  39. mel_feat_conf: Dict = {'n_fft': 1024, 'num_mels': 80, 'sampling_rate': 22050,
  40. 'hop_size': 256, 'win_size': 1024, 'fmin': 0, 'fmax': 8000}):
  41. super().__init__()
  42. self.input_size = input_size
  43. self.output_size = output_size
  44. self.decoder_conf = decoder_conf
  45. self.mel_feat_conf = mel_feat_conf
  46. self.vocab_size = vocab_size
  47. self.output_type = output_type
  48. self.input_frame_rate = input_frame_rate
  49. logging.info(f"input frame rate={self.input_frame_rate}")
  50. self.input_embedding = nn.Embedding(vocab_size, input_size)
  51. self.spk_embed_affine_layer = torch.nn.Linear(spk_embed_dim, output_size)
  52. self.encoder = encoder
  53. self.encoder_proj = torch.nn.Linear(self.encoder.output_size(), output_size)
  54. self.decoder = decoder
  55. self.length_regulator = length_regulator
  56. self.only_mask_loss = only_mask_loss
  57. def forward(
  58. self,
  59. batch: dict,
  60. device: torch.device,
  61. ) -> Dict[str, Optional[torch.Tensor]]:
  62. token = batch['speech_token'].to(device)
  63. token_len = batch['speech_token_len'].to(device)
  64. feat = batch['speech_feat'].to(device)
  65. feat_len = batch['speech_feat_len'].to(device)
  66. embedding = batch['embedding'].to(device)
  67. # xvec projection
  68. embedding = F.normalize(embedding, dim=1)
  69. embedding = self.spk_embed_affine_layer(embedding)
  70. # concat text and prompt_text
  71. mask = (~make_pad_mask(token_len)).float().unsqueeze(-1).to(device)
  72. token = self.input_embedding(torch.clamp(token, min=0)) * mask
  73. # text encode
  74. h, h_lengths = self.encoder(token, token_len)
  75. h = self.encoder_proj(h)
  76. h, h_lengths = self.length_regulator(h, feat_len)
  77. # get conditions
  78. conds = torch.zeros(feat.shape, device=token.device)
  79. for i, j in enumerate(feat_len):
  80. if random.random() < 0.5:
  81. continue
  82. index = random.randint(0, int(0.3 * j))
  83. conds[i, :index] = feat[i, :index]
  84. conds = conds.transpose(1, 2)
  85. mask = (~make_pad_mask(feat_len)).to(h)
  86. feat = F.interpolate(feat.unsqueeze(dim=1), size=h.shape[1:], mode="nearest").squeeze(dim=1)
  87. loss, _ = self.decoder.compute_loss(
  88. feat.transpose(1, 2).contiguous(),
  89. mask.unsqueeze(1),
  90. h.transpose(1, 2).contiguous(),
  91. embedding,
  92. cond=conds
  93. )
  94. return {'loss': loss}
  95. @torch.inference_mode()
  96. def inference(self,
  97. token,
  98. token_len,
  99. prompt_token,
  100. prompt_token_len,
  101. prompt_feat,
  102. prompt_feat_len,
  103. embedding,
  104. required_cache_size=0,
  105. flow_cache=None):
  106. assert token.shape[0] == 1
  107. # xvec projection
  108. embedding = F.normalize(embedding, dim=1)
  109. embedding = self.spk_embed_affine_layer(embedding)
  110. # concat text and prompt_text
  111. token_len1, token_len2 = prompt_token.shape[1], token.shape[1]
  112. token, token_len = torch.concat([prompt_token, token], dim=1), prompt_token_len + token_len
  113. mask = (~make_pad_mask(token_len)).unsqueeze(-1).to(embedding)
  114. token = self.input_embedding(torch.clamp(token, min=0)) * mask
  115. # text encode
  116. h, h_lengths = self.encoder(token, token_len)
  117. h = self.encoder_proj(h)
  118. mel_len1, mel_len2 = prompt_feat.shape[1], int(token_len2 / self.input_frame_rate * 22050 / 256)
  119. h, h_lengths = self.length_regulator.inference(h[:, :token_len1], h[:, token_len1:], mel_len1, mel_len2, self.input_frame_rate)
  120. # get conditions
  121. conds = torch.zeros([1, mel_len1 + mel_len2, self.output_size], device=token.device)
  122. conds[:, :mel_len1] = prompt_feat
  123. conds = conds.transpose(1, 2)
  124. mask = (~make_pad_mask(torch.tensor([mel_len1 + mel_len2]))).to(h)
  125. feat, flow_cache = self.decoder(
  126. mu=h.transpose(1, 2).contiguous(),
  127. mask=mask.unsqueeze(1),
  128. spks=embedding,
  129. cond=conds,
  130. n_timesteps=10,
  131. required_cache_size=required_cache_size,
  132. flow_cache=flow_cache
  133. )
  134. feat = feat[:, :, mel_len1:]
  135. assert feat.shape[2] == mel_len2
  136. return feat, flow_cache