Skip to content

Commit d1c354e

Browse files
committed
add huggingface to pretrained
1 parent b048a2d commit d1c354e

File tree

12 files changed

+2340
-0
lines changed

12 files changed

+2340
-0
lines changed
Lines changed: 6 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,6 @@
1+
FROM verlai/verl:app-verl0.4-vllm0.8.5-mcore0.12.2-te2.2
2+
COPY requirements-cosyvoice.txt /myworkspace/requirements.txt
3+
RUN pip install -r /myworkspace/requirements.txt
4+
RUN pip install -U nvidia-pytriton
5+
RUN git clone https://github.com/yuekaizhang/verl.git /myworkspace/verl -b thread && cd /myworkspace/verl && pip install --no-deps -e .
6+
RUN git clone https://github.com/yuekaizhang/PytritonSenseVoice.git /myworkspace/PytritonSenseVoice && cd /myworkspace/PytritonSenseVoice && pip install -e .

examples/grpo/cosyvoice2/README.md

Lines changed: 121 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,121 @@
1+
# CosyVoice2 LLM Reinforcement Learning Recipe
2+
3+
This recipe demonstrates how to fine-tune the **CosyVoice2** large language model with reinforcement learning algorithms—specifically **GRPO**—using the [veRL](https://github.com/volcengine/verl) framework. Our experiments show that applying GRPO reduces the character error rate (CER) on the CosyVoice3 `zero_shot_zh` set from 4.08 % to 3.36 %.
4+
5+
## Table of Contents
6+
7+
- [Environment Setup](#environment-setup)
8+
- [Data Preparation](#data-preparation)
9+
- [Reward Function & ASR Server](#reward-function--asr-server)
10+
- [Training](#training)
11+
- [Evaluation](#evaluation)
12+
- [Export Model](#export-model)
13+
- [Results](#results)
14+
- [Acknowledgement](#acknowledgement)
15+
16+
## Environment Setup
17+
We recommend using the pre-built Docker image below. Alternatively, you can manually install the dependencies following the Dockerfile.
18+
```bash
19+
docker pull soar97/verl:app-verl0.4-vllm0.8.5-mcore0.12.2-te2.2
20+
```
21+
22+
## Data Preparation
23+
24+
`prepare_data.py` expects a JSON/JSONL file with at least the following schema:
25+
26+
```jsonc
27+
{
28+
"text": "An example sentence to be synthesized."
29+
}
30+
```
31+
You can download the JSONL files from the metadata directory of the [SparkAudio/voxbox](https://huggingface.co/datasets/SparkAudio/voxbox/tree/main/metadata) dataset on Hugging Face.
32+
33+
Stage `0` converts raw JSONL files into the parquet format expected by veRL:
34+
35+
```bash
36+
bash run.sh 0 0
37+
```
38+
Create two JSONL files—`train.jsonl` and `test.jsonl`.
39+
The script will then generate two Parquet files:
40+
41+
```
42+
data/parquet_tiny/train.parquet
43+
data/parquet_tiny/test.parquet
44+
```
45+
46+
Each sample is automatically wrapped into a cosyvoice2-style prompt so that the LLM learns to output CosyVoice2 speech tokens.
47+
48+
49+
## Reward Function & ASR Server
50+
51+
To compute rewards we run a lightweight server that:
52+
53+
1. Converts generated speech tokens back to a 16 kHz waveform with the **CosyVoice2** pretrained U-Net model.
54+
2. Transcribes the waveform with **SenseVoice** ASR.
55+
3. Calculates the pinyin-level error rate relative to the ground-truth text and maps it to a score in the range \[0-1\].
56+
57+
Start the server (stage `1`) in a dedicated terminal or on a separate GPU:
58+
59+
```bash
60+
bash run.sh 1 1
61+
# Triton server listens on ports 8000/8001/8002
62+
```
63+
64+
The custom reward implementation lives in [`reward_tts.py`](./reward_tts.py) and calls the server to obtain the reward score.
65+
66+
## Training
67+
68+
Run stage `2` to start GRPO training:
69+
70+
```bash
71+
bash run.sh 2 2
72+
```
73+
74+
Key CLI arguments passed to `verl.trainer.main_ppo`:
75+
76+
* `algorithm.adv_estimator=grpo` – use GRPO instead of PPO.
77+
* `data.train_files=data/parquet_aishell3/train.parquet` and `data.val_files=data/parquet_aishell3/test.parquet`
78+
* `custom_reward_function.path=reward_tts.py` – custom reward function described above.
79+
80+
Adjust `CUDA_VISIBLE_DEVICES`, batch sizes, and other hyperparameters to match your hardware.
81+
82+
## Evaluation
83+
84+
After training completes, collect the sharded FSDP weights and export a Hugging Face-style checkpoint (stage `3`):
85+
86+
```bash
87+
bash run.sh 3 3 # merges weights into $llm_path/merged_hf_model
88+
```
89+
90+
You can then evaluate the model on the CosyVoice3 zero-shot Chinese test set (stage `4`):
91+
92+
```bash
93+
bash run.sh 4 4
94+
```
95+
96+
This command launches distributed inference via `infer_dataset.py` and computes WER with `scripts/compute_wer.sh`.
97+
98+
> [!TIP]
99+
> The script also supports the Seed-TTS test set by setting `dataset=test_zh`.
100+
101+
## Export Model
102+
103+
To use the RL-trained model with the official CosyVoice repository:
104+
105+
```bash
106+
bash run.sh 5 5
107+
```
108+
109+
The script converts the Hugging Face checkpoint back into the format expected by the CosyVoice repository.
110+
111+
## Results
112+
113+
| Model | Seed-TTS `test_zh` CER | CosyVoice3 `zero_shot_zh` CER | Comment |
114+
|-------|------------------------|------------------------------|---------|
115+
| CosyVoice2 LLM (official) | 1.45 % | 4.08 % | See the [paper](https://arxiv.org/abs/2412.10117) |
116+
| CosyVoice2 LLM + GRPO | 1.37 % | **3.36 %** | See the [decoding results](yuekai/official-cosyvoice-llm-grpo-aishell3) |
117+
118+
## Acknowledgement
119+
120+
This work was inspired by the implementation in [ch-tts-llasa-rl-grpo](https://github.com/channel-io/ch-tts-llasa-rl-grpo).
121+
Lines changed: 70 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,70 @@
1+
2+
# SPDX-FileCopyrightText: Copyright (c) 2025, NVIDIA CORPORATION. All rights reserved.
3+
# SPDX-License-Identifier: Apache-2.0
4+
#
5+
# Licensed under the Apache License, Version 2.0 (the "License");
6+
# you may not use this file except in compliance with the License.
7+
# You may obtain a copy of the License at
8+
#
9+
# http://www.apache.org/licenses/LICENSE-2.0
10+
#
11+
# Unless required by applicable law or agreed to in writing, software
12+
# distributed under the License is distributed on an "AS IS" BASIS,
13+
# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
14+
# See the License for the specific language governing permissions and
15+
# limitations under the License.
16+
"""
17+
python3 hf2pretrained.py --hf-cosyvoice2-llm-path /workspace/rl-exp/checkpoint-400 --output-path /workspace/CosyVoice2-0.5B/llm-new.pt
18+
"""
19+
from argparse import ArgumentParser
20+
import torch
21+
from safetensors import safe_open
22+
from transformers import AutoTokenizer
23+
24+
def get_args():
25+
parser = ArgumentParser()
26+
27+
parser.add_argument(
28+
"--hf-cosyvoice2-llm-path",
29+
type=str,
30+
default=None,
31+
help="The RL trained CosyVoice2 model path in HuggingFace format",
32+
)
33+
parser.add_argument(
34+
"--output-path",
35+
type=str,
36+
default="./llm.pt",
37+
help="The path to save the llm.pt",
38+
)
39+
args = parser.parse_args()
40+
return args
41+
42+
if __name__ == "__main__":
43+
args = get_args()
44+
45+
tokenizer = AutoTokenizer.from_pretrained(args.hf_cosyvoice2_llm_path)
46+
speech_start_idx = tokenizer.convert_tokens_to_ids("<|s_0|>")
47+
cosyvoice2_token_size = 6561 + 3
48+
llm_embedding_vocab_size = 2
49+
50+
hf_tensors = {}
51+
with safe_open(f"{args.hf_cosyvoice2_llm_path}/model.safetensors", framework="pt", device="cpu") as f:
52+
for k in f.keys():
53+
if k.startswith("lm_head.bias"):
54+
# RL trained model disable bias for lm_head
55+
continue
56+
new_k = "llm.model." + k
57+
hf_tensors[new_k] = f.get_tensor(k)
58+
if k.startswith("lm_head"):
59+
hf_tensors["llm_decoder.weight"] = f.get_tensor(k)[speech_start_idx:speech_start_idx + cosyvoice2_token_size]
60+
hf_tensors["llm_decoder.bias"] = torch.zeros_like(hf_tensors["llm_decoder.weight"][:, 0])
61+
if k.startswith("model.embed_tokens"):
62+
hf_tensors["speech_embedding.weight"] = f.get_tensor(k)[speech_start_idx:speech_start_idx + cosyvoice2_token_size]
63+
hf_tensors["llm_embedding.weight"] = f.get_tensor(k)[speech_start_idx + cosyvoice2_token_size:speech_start_idx + cosyvoice2_token_size + llm_embedding_vocab_size]
64+
65+
# use tie_word_embeddings=True
66+
hf_tensors["llm.model.model.embed_tokens.weight"] = hf_tensors["llm.model.model.embed_tokens.weight"][:151936]
67+
hf_tensors["llm.model.lm_head.weight"] = hf_tensors["llm.model.model.embed_tokens.weight"]
68+
69+
torch.save(hf_tensors, args.output_path)
70+

0 commit comments

Comments
 (0)