Skip to content

Commit

Permalink
accumulated update:
Browse files Browse the repository at this point in the history
many dsp operator
accumulate gradient
update sequence length after trimming
filtering in waveglow
metrics for Bayesian model
other loss functions for contrastive loss
...
  • Loading branch information
TonyWangX committed Aug 31, 2022
1 parent 30b3d84 commit c6ec86b
Show file tree
Hide file tree
Showing 32 changed files with 2,310 additions and 130 deletions.
4 changes: 2 additions & 2 deletions README.md
Original file line number Diff line number Diff line change
Expand Up @@ -257,8 +257,8 @@ More instructions can be found in the Jupyter notebook [./tutorials/c01_data_for

Name | Function
------------ | -------------
./core_scripts | scripts to manage the training process, data io, and so on
./core_modules | finished pytorch modules
./core_scripts | scripts (Numpy or Pytorch code) to manage the training process, data io, etc.
./core_modules | finalized pytorch modules
./sandbox | new functions and modules to be test
./project | project directories, and each folder correspond to one model for one dataset
./project/\*/\*/main.py | script to load data and run training and inference
Expand Down
42 changes: 35 additions & 7 deletions core_scripts/config_parse/arg_parse.py
Original file line number Diff line number Diff line change
Expand Up @@ -43,6 +43,10 @@ def f_args_parsed(argument_input = None):
# Training settings
mes = 'batch size for training/inference (default: 1)'
parser.add_argument('--batch-size', type=int, default=1, help=mes)


mes = 'number of mini-batches to accumulate (default: 1)'
parser.add_argument('--size-accumulate-grad', type=int, default=1, help=mes)

mes = 'number of epochs to train (default: 50)'
parser.add_argument('--epochs', type=int, default=50, help=mes)
Expand All @@ -68,6 +72,19 @@ def f_args_parsed(argument_input = None):
mes += ' Training stopped after --no-best-epochs.'
parser.add_argument('--lr-decay-factor', type=float, default=-1.0, help=mes)

mes = 'lr scheduler: 0: ReduceLROnPlateau (default); 1: StepLR; '
mes += 'this option is set on only when --lr-decay-factor > 0. '
mes += 'Please check core_scripts/op_manager/lr_scheduler.py '
mes += 'for detailed hyper config for each type of lr scheduler'
parser.add_argument('--lr-scheduler-type', type=int, default=0, help=mes)

mes = 'lr patience: patience for torch_optim_steplr.ReduceLROnPlateau '
mes += 'this option is used only when --lr-scheduler-type == 0. '
parser.add_argument('--lr-patience', type=int, default=5, help=mes)

mes = 'lr step size: step size for torch.optim.lr_scheduler.StepLR'
mes += 'this option is used only when --lr-scheduler-type == 1. '
parser.add_argument('--lr-steplr-size', type=int, default=5, help=mes)

mes = 'L2 penalty on weight (default: not use). '
mes += 'It corresponds to the weight_decay option in Adam'
Expand All @@ -77,13 +94,7 @@ def f_args_parsed(argument_input = None):
mes += 'default (-1, not use)'
parser.add_argument('--grad-clip-norm', type=float, default=-1.0,
help=mes)

mes = 'lr scheduler: 0: ReduceLROnPlateau (default); 1: StepLR; '
mes += 'this option is set on only when --lr-decay-factor > 0. '
mes += 'Please check core_scripts/op_manager/lr_scheduler.py '
mes += 'for detailed hyper config for each type of lr scheduler'
parser.add_argument('--lr-scheduler-type', type=int, default=0, help=mes)


parser.add_argument('--no-cuda', action='store_true', default=False,
help='disables CUDA training')

Expand Down Expand Up @@ -134,6 +145,10 @@ def f_args_parsed(argument_input = None):
mes = "Ignore existing cache file dic"
parser.add_argument('--ignore-cached-file-infor',
action='store_true', default=False, help=mes)

mes = "External directory to store cache file dic"
parser.add_argument('--path-cache-file', type=str, default="", help=mes)

######
# options to save model / checkpoint
parser.add_argument('--save-model-dir', type=str, \
Expand Down Expand Up @@ -223,6 +238,11 @@ def f_args_parsed(argument_input = None):
mes = 'path to save generated data (default: ./output)'
parser.add_argument('--output-dir', type=str, default="./output", \
help=mes)

# options to output
mes = 'prefix added to file name (default: no string)'
parser.add_argument('--output-filename-prefix', type=str, default="", \
help=mes)

mes = 'truncate input data sequences so that the max length < N.'
mes += ' (default: -1, not do truncating at all)'
Expand Down Expand Up @@ -299,6 +319,14 @@ def f_args_parsed(argument_input = None):
default=0, help=mes)


mes = 'update data length in internal buffer if data length is changed '
mes += 'by augmentation method. This is useful, for example, when using '
mes += '--sampler block_shuffle_by_length --opt-wav-silence-handler 3 '
mes += 'or using other data augmentation method changes data length.'
parser.add_argument('--force-update-seq-length', action='store_true', \
default=False, help=mes)


#
# done
if argument_input is not None:
Expand Down
1 change: 1 addition & 0 deletions core_scripts/data_io/customize_collate_fn.py
Original file line number Diff line number Diff line change
Expand Up @@ -114,6 +114,7 @@ def customize_collate(batch):
numel = max([x.numel() for x in batch_new]) * len(batch_new)
storage = elem.storage()._new_shared(numel)
out = elem.new(storage)
#print(batch_new.shape[0], batch_new.shape[1])
return torch.stack(batch_new, 0, out=out)

elif elem_type.__module__ == 'numpy' and elem_type.__name__ != 'str_' \
Expand Down
66 changes: 63 additions & 3 deletions core_scripts/data_io/customize_dataset.py
Original file line number Diff line number Diff line change
Expand Up @@ -22,6 +22,7 @@
import core_scripts.data_io.customize_collate_fn as nii_collate_fn
import core_scripts.data_io.customize_sampler as nii_sampler_fn
import core_scripts.data_io.conf as nii_dconf
import core_scripts.data_io.seq_info as nii_seqinfo

__author__ = "Xin Wang"
__email__ = "[email protected]"
Expand All @@ -35,7 +36,7 @@
class merge_loader():
""" merge_loader
Data loader for customized data with m_concate_set = None
By defauly, draw equal number of samples from each subset
By default, draw equal number of samples from each subset
__iter__():
__next__(): load data and merge into minibatch
Expand Down Expand Up @@ -162,6 +163,31 @@ def f_get_seq_len_list(self):
tmp += sub_dataset.f_get_seq_len_list()
return tmp

def f_get_updated_seq_len_for_sampler_list(self):
""" Similar to f_get_seq_len_list
but it returns the updated data sequence length only for
length-based shuffling in sampler
"""
tmp = []
for sub_dataset in self.datasets:
tmp += sub_dataset.f_get_updated_seq_len_for_sampler_list()
return tmp

def f_update_seq_len_for_sampler_list(self, data_info):
"""
"""
for one_info in data_info:
data_idx = nii_seqinfo.parse_idx(one_info)
data_len = nii_seqinfo.parse_length(one_info)
for idx_u, idx_d, subset in \
zip(self.len_top, self.len_bot, self.datasets):
if data_idx < idx_u:
subset.f_update_seq_len_for_sampler_list(data_idx, data_len)
break
else:
pass
return

def f_manage_data(self, lst_data_idx, opt):
""" f_manage_data(self, lst_data_idx, opt)
"""
Expand Down Expand Up @@ -420,13 +446,14 @@ def print_info(self):
dset.print_info()
return

def putitem(self, output_data, save_dir, data_infor_str):
def putitem(self, output_data, save_dir, filename_prefix, data_infor_str):
""" Decompose the output_data from network into
separate files
"""
# Since all datasets have similar configuration on feat dim,
# use anyone is OK
self.m_datasets[0].putitem(output_data, save_dir, data_infor_str)
self.m_datasets[0].putitem(output_data, save_dir, filename_prefix,
data_infor_str)

def get_in_dim(self):
""" Return the dimension of input features
Expand Down Expand Up @@ -457,6 +484,39 @@ def get_seq_list(self):
tmp += dataset.get_seq_list()
return tmp

def update_seq_len_in_sampler_sub(self, data_info):
"""
"""
# assume data_info logs the new data length that can be used for
# sampler shuffle_by_length
if self.way_to_merge == 'concatenate':
self.m_concate_set.f_update_seq_len_for_sampler_list(data_info)
else:
print("Not implemented")
sys.exit(1)
return

def update_seq_len_in_sampler(self):
""" update_seq_len()
Update sequence length if sequence length has been changed
(for example, during silence trim process)
This is necessary when using shuffle_by_seq_length sampler
and the sequences were trimmed in data augmentation function.
"""
# call each subdataset and update the sequence length
for idx, _ in enumerate(self.m_datasets):
self.m_datasets[idx].update_seq_len_in_sampler()

# update loader of this database
if self.way_to_merge == 'concatenate':
if self.m_params['sampler'] == nii_sampler_fn.g_str_sampler_bsbl \
and hasattr(self.m_loader.sampler, 'update_seq_length'):
self.m_loader.sampler.update_seq_length(
self.m_concate_set.f_get_updated_seq_len_for_sampler_list())
return

def manage_data(self, lst_data_idx, opt):
""" manage_data(data_index_list, opt)
Expand Down
12 changes: 12 additions & 0 deletions core_scripts/data_io/customize_sampler.py
Original file line number Diff line number Diff line change
Expand Up @@ -76,6 +76,18 @@ def __len__(self):
https://pytorch.org/docs/stable/data.html#torch.utils.data.Sampler
"""
return len(self.m_idx)

def update_seq_length(self, buf_dataseq_length):
"""Update sequence length if necessary
This will resort the sequences based on updated sequence length
"""
if len(buf_dataseq_length) == len(self.m_idx):
self.m_idx = np.argsort(buf_dataseq_length)
else:
print("Incompatible sequence length input: updata_seq_length")
sys.exit(1)
return

if __name__ == "__main__":
print("Definition of customized_sampler")
Loading

0 comments on commit c6ec86b

Please sign in to comment.