6
6
import wandb
7
7
from coati .distributed .consumer import BaseConsumer
8
8
from coati .distributed .loss import PolicyLoss
9
- from coati .distributed .utils import memory_efficient_logprob
9
+ from coati .distributed .utils import entropy_from_logits , memory_efficient_logprob
10
10
from coati .trainer .utils import all_reduce_mean , all_reduce_sum
11
11
from transformers import AutoModelForCausalLM , AutoTokenizer
12
12
@@ -75,6 +75,7 @@ def __init__(
75
75
self .optimizer = HybridAdam (self .policy_model .parameters (), lr = grpo_config .get ("lr" , 1e-6 ))
76
76
self .accum_loss = torch .zeros (1 , device = self .device )
77
77
self .accum_kl = torch .zeros (1 , device = self .device )
78
+ self .accum_entropy = torch .zeros (1 , device = self .device )
78
79
self .accum_advantages = torch .zeros (1 , device = self .device )
79
80
self .raw_train_batch_reward = []
80
81
self .raw_train_batch_format_acc = []
@@ -244,6 +245,7 @@ def step(self, step_idx: int, pbar: Any, **kwargs) -> Optional[float]:
244
245
else self .booster .no_sync (self .policy_model , self .optimizer )
245
246
)
246
247
with ctx :
248
+ mini_batch_entropies = []
247
249
for forward_micro_batch_start in range (0 , data ["input_ids" ].size (0 ), train_microbatch_size ):
248
250
input_ids_forward_micro_batch = data ["input_ids" ][
249
251
forward_micro_batch_start : forward_micro_batch_start + train_microbatch_size
@@ -310,9 +312,11 @@ def step(self, step_idx: int, pbar: Any, **kwargs) -> Optional[float]:
310
312
data_policy_forward ["reference_action_log_probs" ] = reference_action_log_probs
311
313
312
314
kl = []
315
+ policy_model_logits = torch .empty_like (input_ids_forward_micro_batch , device = self .device )
313
316
314
317
def _criterion (outputs , inputs ):
315
318
action_logits = outputs .logits
319
+ policy_model_logits .copy_ (action_logits )
316
320
action_log_probs = memory_efficient_logprob (
317
321
action_logits / self .generate_config ["temperature" ],
318
322
inputs ["input_ids" ],
@@ -359,6 +363,20 @@ def _criterion(outputs, inputs):
359
363
kl = all_reduce_mean (torch .mean (torch .stack (kl )).to (loss .device ), self .plugin ).data
360
364
mean_kl .append (kl )
361
365
mean_loss .append (all_reduce_mean (loss , self .plugin ).data )
366
+ mini_batch_entropies .append (
367
+ all_reduce_mean (
368
+ (
369
+ (
370
+ (
371
+ entropy_from_logits (policy_model_logits [:, - num_action :])
372
+ * action_mask_forward_micro_batch
373
+ ).sum (- 1 )
374
+ )
375
+ / action_mask_forward_micro_batch .sum (- 1 )
376
+ ).detach (),
377
+ self .plugin ,
378
+ )
379
+ )
362
380
else :
363
381
policy_model_logits = self .policy_model (
364
382
input_ids = input_ids_forward_micro_batch ,
@@ -412,6 +430,20 @@ def _criterion(outputs, inputs):
412
430
kl = all_reduce_mean (kl .mean (), self .plugin )
413
431
mean_kl .append (kl .data )
414
432
mean_loss .append (loss .data )
433
+ mini_batch_entropies .append (
434
+ all_reduce_mean (
435
+ (
436
+ (
437
+ (
438
+ entropy_from_logits (policy_model_logits [:, - num_action :])
439
+ * action_mask_forward_micro_batch
440
+ ).sum (- 1 )
441
+ )
442
+ / action_mask_forward_micro_batch .sum (- 1 )
443
+ ).detach (),
444
+ self .plugin ,
445
+ )
446
+ )
415
447
if not self .plugin .pp_size > 1 or (
416
448
self .plugin .pp_size > 1
417
449
and self .booster .plugin .stage_manager .is_last_stage ()
@@ -423,7 +455,9 @@ def _criterion(outputs, inputs):
423
455
ans_acc = all_reduce_mean (ans_acc .mean (), self .plugin )
424
456
advantages = all_reduce_mean (advantages .mean (), self .plugin )
425
457
response_length = all_reduce_mean (response_length .mean (), self .plugin )
458
+ entropy = torch .cat (mini_batch_entropies , dim = 0 ).mean ()
426
459
self .accum_loss .add_ (sum (mean_loss ) / len (mean_loss ))
460
+ self .accum_entropy .add_ (entropy .data )
427
461
if self .policy_loss_fn .beta > 0 :
428
462
self .accum_kl .add_ (sum (mean_kl ) / len (mean_kl ))
429
463
self .accum_advantages .add_ (advantages .data )
@@ -464,6 +498,7 @@ def _criterion(outputs, inputs):
464
498
f"Response Length: { raw_batch_response_len_mean :.4f} " ,
465
499
f"Sample_utilization: { sample_utilization :.4f} " ,
466
500
f"Overlength samples ratio: { overlength_samples_ratio :.4f} " ,
501
+ f"Entropy: { self .accum_entropy .item () / self .accum_count :.4f} " ,
467
502
] + ([f"KL: { self .accum_kl .item () / self .accum_count :.4f} " ] if self .policy_loss_fn .beta > 0 else [])
468
503
print ("\n " .join (to_log_msg ))
469
504
metrics = {
@@ -475,6 +510,7 @@ def _criterion(outputs, inputs):
475
510
"train/advantages" : self .accum_advantages .item () / self .accum_count ,
476
511
"train/learning_rate" : self .lr_scheduler .get_last_lr ()[0 ],
477
512
"train/sample_utilization" : sample_utilization ,
513
+ "train/entropy" : self .accum_entropy .item () / self .accum_count ,
478
514
"train/overlength_samples_ratio" : overlength_samples_ratio ,
479
515
"rollout/temperature" : data ["temperature" ].cpu ().numpy ()[0 ][0 ],
480
516
}
@@ -484,6 +520,7 @@ def _criterion(outputs, inputs):
484
520
self .wandb_run .log (metrics )
485
521
self .accum_loss .zero_ ()
486
522
self .accum_kl .zero_ ()
523
+ self .accum_entropy .zero_ ()
487
524
self .accum_advantages .zero_ ()
488
525
self .accum_count = 0
489
526
return loss_scalar
0 commit comments