Skip to content

Commit fb61428

Browse files
authored
Add tolerable_loss to TuningConfig (#1579)
Signed-off-by: Kaihui-intel <[email protected]>
1 parent 7bf89eb commit fb61428

File tree

5 files changed

+205
-5
lines changed

5 files changed

+205
-5
lines changed

neural_compressor/common/base_tuning.py

+66-3
Original file line numberDiff line numberDiff line change
@@ -213,14 +213,57 @@ class TuningConfig:
213213
config_set: quantization configs. Default value is empty.
214214
timeout: Tuning timeout (seconds). Default value is 0 which means early stop.
215215
max_trials: Max tuning times. Default value is 100. Combine with timeout field to decide when to exit.
216+
tolerable_loss: This float indicates how much metric loss we can accept. \
217+
The metric loss is relative, it can be both positive and negative. Default is 0.01.
218+
219+
Examples:
220+
from neural_compressor import TuningConfig
221+
tune_config = TuningConfig(
222+
config_set=[config1, config2, ...],
223+
max_trials=3,
224+
tolerable_loss=0.01
225+
)
226+
227+
# Case 1: Tolerable Loss
228+
fp32_baseline = 100
229+
config1_metric, config2_metric, ... = 98, 99, ...
230+
231+
# Tuning result of case 1:
232+
# The best tuning config is config2, because config2_metric >= fp32_baseline * (1 - tolerable_loss)
233+
234+
# Case 2: Maximum Trials
235+
fp32_baseline = 100
236+
config1_metric, config2_metric, config3_metric, ... = 98, 98, 97, ...
237+
238+
# Tuning result of case 2:
239+
# The best tuning config is config2, because of the following:
240+
# 1. Not achieving the set goal. (config_metric < fp32_baseline * (1 - tolerable_loss))
241+
# 2. Reached maximum tuning times.
242+
243+
# Case 3: Timeout
244+
tune_config = TuningConfig(
245+
config_set=[config1, config2, ...],
246+
timeout=10, # seconds
247+
max_trials=3,
248+
tolerable_loss=0.01
249+
)
250+
config1_tuning_time, config2_tuning_time, config3_tuning_time, ... = 4, 5, 6, ... # seconds
251+
fp32_baseline = 100
252+
config1_metric, config2_metric, config3_metric, ... = 98, 98, 97, ...
253+
254+
# Tuning result of case 3:
255+
# The best tuning config is config2, due to timeout, the third trial was forced to exit.
216256
"""
217257

218-
def __init__(self, config_set=None, timeout=0, max_trials=100, sampler: Sampler = None) -> None:
258+
def __init__(
259+
self, config_set=None, timeout=0, max_trials=100, sampler: Sampler = None, tolerable_loss=0.01
260+
) -> None:
219261
"""Init a TuneCriterion object."""
220262
self.config_set = config_set
221263
self.timeout = timeout
222264
self.max_trials = max_trials
223265
self.sampler = sampler
266+
self.tolerable_loss = tolerable_loss
224267

225268

226269
class _TrialRecord:
@@ -242,12 +285,17 @@ def __init__(self, tuning_config: TuningConfig) -> None:
242285
self.tuning_config = tuning_config
243286
self.trial_cnt = 0
244287
self.tuning_history: List[_TrialRecord] = []
288+
self.baseline = None
245289

246290
def add_trial_result(self, trial_index: int, trial_result: Union[int, float], quant_config: BaseConfig) -> None:
247291
self.trial_cnt += 1
248292
trial_record = _TrialRecord(trial_index, trial_result, quant_config)
249293
self.tuning_history.append(trial_record)
250294

295+
def set_baseline(self, baseline: float):
296+
self.baseline = baseline
297+
logger.info(f"Fp32 baseline is {self.baseline}")
298+
251299
def get_number_of_trials(self):
252300
return len(self.tuning_history)
253301

@@ -260,8 +308,23 @@ def get_best_quant_config(self) -> BaseConfig:
260308
return sorted_trials_records[0].quant_config
261309

262310
def need_stop(self) -> bool:
263-
# TODO Support more stop criteria in the next PR, such as `reach accuracy goal`, `timeout`, and so on.
264-
return self.trial_cnt >= self.tuning_config.max_trials
311+
"""Check if need to stop tuning. Either accuracy goal is met, max trials is reached or timeout is reached.
312+
313+
Returns:
314+
bool: True if need to stop, otherwise False.
315+
"""
316+
317+
# TODO: Support more stop criteria in the next PR, such as `timeout`, and so on.
318+
# reach max trials
319+
reach_max_trials = self.trial_cnt >= self.tuning_config.max_trials
320+
# reach accuracy goal
321+
meet_accuracy_goal = (
322+
False
323+
if self.baseline is None
324+
else self.tuning_history[-1].trial_result >= (self.baseline * (1 - self.tuning_config.tolerable_loss))
325+
)
326+
# [-1] is the last element representing the latest trail record.
327+
return reach_max_trials or meet_accuracy_goal
265328

266329

267330
def init_tuning(tuning_config: TuningConfig) -> Tuple[ConfigLoader, TuningLogger, TuningMonitor]:

neural_compressor/torch/quantization/autotune.py

+4-1
Original file line numberDiff line numberDiff line change
@@ -46,6 +46,8 @@ def autotune(
4646
evaluator.set_eval_fn_registry(eval_fns)
4747
evaluator.self_check()
4848
config_loader, tuning_logger, tuning_monitor = init_tuning(tuning_config=tune_config)
49+
baseline: float = evaluator.evaluate(model)
50+
tuning_monitor.set_baseline(baseline)
4951
tuning_logger.tuning_start()
5052
for trial_index, quant_config in enumerate(config_loader):
5153
tuning_logger.trial_start(trial_index=trial_index)
@@ -58,11 +60,12 @@ def autotune(
5860
eval_result: float = evaluator.evaluate(q_model)
5961
tuning_logger.evaluation_end()
6062
tuning_monitor.add_trial_result(trial_index, eval_result, quant_config)
63+
tuning_logger.trial_end(trial_index)
6164
if tuning_monitor.need_stop():
6265
best_quant_config: BaseConfig = tuning_monitor.get_best_quant_config()
6366
# !!! Make sure to use deepcopy only when inplace is set to `True`.
6467
quantize(deepcopy(model), quant_config=best_quant_config, run_fn=run_fn, run_args=run_args, inplace=True)
6568
best_quant_model = model # quantize model inplace
66-
tuning_logger.trial_end(trial_index)
69+
break
6770
tuning_logger.tuning_end()
6871
return best_quant_model

test/3x/common/test_common.py

+38-1
Original file line numberDiff line numberDiff line change
@@ -42,7 +42,13 @@
4242

4343
from typing import Any, Callable, List, Optional, Tuple, Union
4444

45-
from neural_compressor.common.base_config import BaseConfig, get_all_config_set_from_config_registry, register_config
45+
from neural_compressor.common.base_config import (
46+
BaseConfig,
47+
ComposableConfig,
48+
get_all_config_set_from_config_registry,
49+
register_config,
50+
)
51+
from neural_compressor.common.base_tuning import ConfigLoader, Sampler
4652
from neural_compressor.common.utils import DEFAULT_WHITE_LIST, OP_NAME_OR_MODULE_TYPE
4753

4854
PRIORITY_FAKE_ALGO = 100
@@ -137,5 +143,36 @@ def test_api(self):
137143
self.assertEqual(config_set[0].weight_bits, DEFAULT_WEIGHT_BITS)
138144

139145

146+
class TestConfigLoader(unittest.TestCase):
147+
def setUp(self):
148+
self.config_set = [get_default_fake_config(), get_default_fake_config()]
149+
self.loader = ConfigLoader(self.config_set, Sampler())
150+
151+
def test_parse_quant_config_single(self):
152+
quant_config = get_default_fake_config()
153+
result = ConfigLoader.parse_quant_config(quant_config)
154+
self.assertEqual(str(result), str(quant_config.expand()))
155+
156+
def test_parse_quant_config_composable(self):
157+
quant_config = get_default_fake_config()
158+
composable_config = ComposableConfig(get_default_fake_config())
159+
composable_config.config_list = [quant_config]
160+
result = ConfigLoader.parse_quant_config(composable_config)
161+
self.assertEqual(str(result), str(quant_config.expand()))
162+
163+
def test_parse_quant_configs(self):
164+
quant_configs = [get_default_fake_config(), get_default_fake_config()]
165+
self.config_set[0].expand = lambda: quant_configs
166+
self.config_set[1].expand = lambda: []
167+
result = self.loader.parse_quant_configs()
168+
self.assertEqual(result, quant_configs)
169+
170+
def test_iteration(self):
171+
quant_configs = [get_default_fake_config(), get_default_fake_config()]
172+
self.loader.parse_quant_configs = lambda: quant_configs
173+
result = list(self.loader)
174+
self.assertEqual(result, quant_configs)
175+
176+
140177
if __name__ == "__main__":
141178
unittest.main()

test/3x/common/test_utility.py

+58
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,58 @@
1+
"""Tests for common components.
2+
3+
!!! Please do not import any framework-specific modules in this file. !!!
4+
* Note, we may need to add some auto check mechanisms to ensure this.
5+
6+
These tests aim to assess the fundamental functionalities of common utils and enhance code coverage.
7+
All tests will be included for each framework CI.
8+
"""
9+
import unittest
10+
11+
from neural_compressor.common import options
12+
from neural_compressor.common.utils import set_random_seed, set_resume_from, set_tensorboard, set_workspace
13+
14+
15+
class TestOptions(unittest.TestCase):
16+
def test_set_random_seed(self):
17+
seed = 12345
18+
set_random_seed(seed)
19+
self.assertEqual(options.random_seed, seed)
20+
21+
# non int type
22+
seed = "12345"
23+
with self.assertRaises(AssertionError):
24+
set_random_seed(seed)
25+
26+
def test_set_workspace(self):
27+
workspace = "/path/to/workspace"
28+
set_workspace(workspace)
29+
self.assertEqual(options.workspace, workspace)
30+
31+
# non String type
32+
workspace = 12345
33+
with self.assertRaises(AssertionError):
34+
set_workspace(workspace)
35+
36+
def test_set_resume_from(self):
37+
resume_from = "/path/to/resume"
38+
set_resume_from(resume_from)
39+
self.assertEqual(options.resume_from, resume_from)
40+
41+
# non String type
42+
resume_from = 12345
43+
with self.assertRaises(AssertionError):
44+
set_resume_from(resume_from)
45+
46+
def test_set_tensorboard(self):
47+
tensorboard = True
48+
set_tensorboard(tensorboard)
49+
self.assertEqual(options.tensorboard, tensorboard)
50+
51+
# non bool type
52+
tensorboard = 123
53+
with self.assertRaises(AssertionError):
54+
set_tensorboard(tensorboard)
55+
56+
57+
if __name__ == "__main__":
58+
unittest.main()

test/3x/torch/test_autotune.py

+39
Original file line numberDiff line numberDiff line change
@@ -200,6 +200,45 @@ def test_autotune_not_eval_func(self):
200200
str(context.exception), "Please ensure that you register at least one evaluation metric for auto-tune."
201201
)
202202

203+
def test_autotune_baseline(self):
204+
logger.info("test_autotune_api")
205+
from neural_compressor.common.base_tuning import evaluator
206+
207+
baseline = [1.0]
208+
209+
# case 1
210+
# Where default tolerable_loss is 0.01, we expect the tuning to end with a "2-trail end" output logged.
211+
acc_res_lst = baseline + [0.9] * 2 + [0.99]
212+
213+
def eval_acc_fn(model):
214+
res = acc_res_lst.pop(0)
215+
return res
216+
217+
custom_tune_config = TuningConfig(config_set=[RTNConfig(bits=[4, 6, 5, 8])], max_trials=6)
218+
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fns=eval_acc_fn)
219+
self.assertIsNotNone(best_model)
220+
221+
# case 2
222+
# Where tolerable_loss is 0.1, we expect the tuning to end with a "0-trail end" output logged.
223+
acc_res_lst = baseline + [0.9] * 2 + [0.99] + [1.01]
224+
custom_tune_config = TuningConfig(config_set=[RTNConfig(bits=[4, 6, 5, 8])], tolerable_loss=0.1)
225+
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fns=eval_acc_fn)
226+
self.assertIsNotNone(best_model)
227+
228+
# case 3
229+
# Where tolerable_loss is -0.01, we expect the tuning to end with a "3-trail end" output logged.
230+
acc_res_lst = baseline + [0.9] * 2 + [0.99] + [1.01]
231+
custom_tune_config = TuningConfig(config_set=[RTNConfig(bits=[4, 6, 5, 8])], tolerable_loss=-0.01)
232+
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fns=eval_acc_fn)
233+
self.assertIsNotNone(best_model)
234+
235+
# case 4
236+
# Where tolerable_loss is 0.01 and accuracy meets the goal, we expect best model is None.
237+
acc_res_lst = baseline + [0.9] * 2 + [0.9] + [0.9]
238+
custom_tune_config = TuningConfig(config_set=[RTNConfig(bits=[4, 6, 5, 8])], tolerable_loss=0.01)
239+
best_model = autotune(model=build_simple_torch_model(), tune_config=custom_tune_config, eval_fns=eval_acc_fn)
240+
self.assertIsNone(best_model)
241+
203242

204243
if __name__ == "__main__":
205244
unittest.main()

0 commit comments

Comments
 (0)