Skip to content

Commit d10c76c

Browse files
xin3hechensuyuepre-commit-ci[bot]
authored
add tuning result table (#2339)
Signed-off-by: He, Xin3 <xin3.he@intel.com> Co-authored-by: chen, suyue <suyue.chen@intel.com> Co-authored-by: pre-commit-ci[bot] <66853113+pre-commit-ci[bot]@users.noreply.github.com>
1 parent cb4662f commit d10c76c

File tree

3 files changed

+56
-6
lines changed

3 files changed

+56
-6
lines changed

neural_compressor/common/base_tuning.py

Lines changed: 48 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -18,7 +18,7 @@
1818
from typing import Any, Callable, Dict, Generator, Iterator, List, Optional, Sized, Tuple, Union
1919

2020
from neural_compressor.common.base_config import BaseConfig
21-
from neural_compressor.common.utils import TuningLogger, logger
21+
from neural_compressor.common.utils import Statistics, TuningLogger, logger
2222

2323
__all__ = [
2424
"Evaluator",
@@ -423,6 +423,47 @@ def add_trial_result(self, trial_index: int, trial_result: Union[int, float], qu
423423
trial_record = _TrialRecord(trial_index, trial_result, quant_config)
424424
self.tuning_history.append(trial_record)
425425

426+
# Print tuning results table
427+
self._print_trial_results_table(trial_index, trial_result)
428+
429+
def _print_trial_results_table(self, trial_index: int, trial_result: Union[int, float]) -> None:
430+
"""Print trial results in a formatted table using Statistics class."""
431+
baseline_val = self.baseline if self.baseline is not None else 0.0
432+
baseline_str = f"{baseline_val:.4f}" if self.baseline is not None else "N/A"
433+
target_threshold_str = (
434+
f"{baseline_val * (1 - self.tuning_config.tolerable_loss):.4f}" if self.baseline is not None else "N/A"
435+
)
436+
437+
# Calculate relative loss if baseline is available
438+
relative_loss_val = 0.0
439+
relative_loss_str = "N/A"
440+
if self.baseline is not None:
441+
relative_loss_val = (baseline_val - trial_result) / baseline_val
442+
relative_loss_str = f"{relative_loss_val*100:.2f}%"
443+
444+
# Get best result so far
445+
best_result = max(record.trial_result for record in self.tuning_history)
446+
447+
# Status indicator with emoji
448+
if self.baseline is not None and trial_result >= (baseline_val * (1 - self.tuning_config.tolerable_loss)):
449+
status = "✅ PASSED"
450+
else:
451+
status = "❌ FAILED"
452+
453+
# Prepare data for Statistics table with combined fields
454+
field_names = ["📊 Metric", "📈 Value"]
455+
output_data = [
456+
["Trial / Progress", f"{len(self.tuning_history)}/{self.tuning_config.max_trials}"],
457+
["Baseline / Target", f"{baseline_str} / {target_threshold_str}"],
458+
["Current / Status", f"{trial_result:.4f} | {status}"],
459+
["Best / Relative Loss", f"{best_result:.4f} / {relative_loss_str}"],
460+
]
461+
462+
# Use Statistics class to print the table
463+
Statistics(
464+
output_data, header=f"🎯 Auto-Tune Trial #{trial_index} Results", field_names=field_names
465+
).print_stat()
466+
426467
def set_baseline(self, baseline: float):
427468
"""Set the baseline value for auto-tune.
428469
@@ -488,4 +529,10 @@ def init_tuning(tuning_config: TuningConfig) -> Tuple[ConfigLoader, TuningLogger
488529
config_loader = ConfigLoader(config_set=tuning_config.config_set, sampler=tuning_config.sampler)
489530
tuning_logger = TuningLogger()
490531
tuning_monitor = TuningMonitor(tuning_config)
532+
533+
# Update max_trials based on actual number of available configurations
534+
actual_config_count = len(config_loader.config_set)
535+
if tuning_config.max_trials > actual_config_count:
536+
tuning_config.max_trials = actual_config_count
537+
491538
return config_loader, tuning_logger, tuning_monitor

test/3x/torch/quantization/weight_only/test_autoround.py

Lines changed: 3 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -370,6 +370,7 @@ def test_scheme(self, scheme):
370370

371371

372372
@pytest.mark.skipif(not ct_installed, reason="The compressed-tensors module is not installed.")
373+
@pytest.mark.skipif(Version(auto_round.__version__) < Version("0.9.0"), reason="target bits is not supported.")
373374
def test_target_bits(self):
374375
fp32_model = AutoModelForCausalLM.from_pretrained(
375376
"facebook/opt-125m",
@@ -401,6 +402,8 @@ def test_target_bits(self):
401402
"model is not quantized correctly, please check."
402403

403404

405+
@pytest.mark.skipif(not ct_installed, reason="The compressed-tensors module is not installed.")
406+
@pytest.mark.skipif(Version(auto_round.__version__) < Version("0.9.0"), reason="target bits is not supported.")
404407
def test_target_bits_autotune(self):
405408
from neural_compressor.torch.quantization import TuningConfig, autotune
406409
baseline = 1

test/3x/torch/test_autotune.py

Lines changed: 5 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -299,11 +299,11 @@ def eval_acc_fn(model):
299299
self.assertIsNotNone(best_model)
300300

301301
# case 4
302-
# Where tolerable_loss is 0.01 and accuracy meets the goal, we expect best model is None.
302+
# Where tolerable_loss is 0.01 and accuracy doesn't meets the goal, best_model is the best model in trails.
303303
acc_res_lst = baseline + [0.9] * 2 + [0.9] + [0.9]
304304
custom_tune_config = TuningConfig(config_set=[RTNConfig(bits=[4, 6, 5, 8])], tolerable_loss=0.01)
305305
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fn=eval_acc_fn)
306-
self.assertIsNone(best_model)
306+
self.assertIsNotNone(best_model)
307307

308308
@reset_tuning_target
309309
def test_rtn_double_quant_config_set(self) -> None:
@@ -335,7 +335,7 @@ def eval_acc_fn(model) -> float:
335335
config_set=get_rtn_double_quant_config_set(), max_trials=10, tolerable_loss=-1
336336
)
337337
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fn=eval_acc_fn)
338-
self.assertIsNone(best_model)
338+
self.assertIsNotNone(best_model)
339339

340340
@patch("neural_compressor.torch.utils.constants.DOUBLE_QUANT_CONFIGS", FAKE_DOUBLE_QUANT_CONFIGS)
341341
def test_rtn_double_quant_config_set3(self) -> None:
@@ -350,7 +350,7 @@ def eval_acc_fn(model) -> float:
350350

351351
custom_tune_config = TuningConfig(config_set=get_rtn_double_quant_config_set(), tolerable_loss=-1)
352352
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fn=eval_acc_fn)
353-
self.assertIsNone(best_model)
353+
self.assertIsNotNone(best_model)
354354

355355
def test_woq_tuning(self):
356356
from neural_compressor.torch.quantization import autotune, get_woq_tuning_config
@@ -374,7 +374,7 @@ def eval_acc_fn(model):
374374
run_args=(dataloader, True), # run_args should be a tuple,
375375
example_inputs=example_inputs,
376376
)
377-
self.assertIsNone(best_model)
377+
self.assertIsNotNone(best_model)
378378

379379
@reset_tuning_target
380380
def test_autotune_mixed_precision_default(self):

0 commit comments

Comments
 (0)